Skip to content

Commit

Permalink
feat: 🎸 add a LTSF baseline SOFTS (added by @superarthurlx)
Browse files Browse the repository at this point in the history
  • Loading branch information
zezhishao committed Nov 27, 2024
1 parent 47e0070 commit c2940f2
Show file tree
Hide file tree
Showing 6 changed files with 307 additions and 10 deletions.
1 change: 1 addition & 0 deletions README.md
Original file line number Diff line number Diff line change
Expand Up @@ -134,6 +134,7 @@ The code links (💻Code) in the table below point to the official implementatio
| 📊Baseline | 📝Title | 📄Paper | 💻Code | 🏛Venue | 🎯Task |
| :------------ | :------------------------------------------------------------------------------------------------------- | :----------------------------------------------------- | :---------------------------------------------------------------------------- | :--------- | :----- |
| SOFTS | SOFTS: Efficient Multivariate Time Series Forecasting with Series-Core Fusion | [Link](https://arxiv.org/pdf/2404.14197) | [Link](https://github.com/Secilia-Cxy/SOFTS) | NeurIPS'24 | LTSF |
| CATS | Are Self-Attentions Effective for Time Series Forecasting? | [Link](https://arxiv.org/pdf/2405.16877) | [Link](https://github.com/dongbeank/CATS) | NeurIPS'24 | LTSF |
| Sumba | Structured Matrix Basis for Multivariate Time Series Forecasting with Interpretable Dynamics | [Link](https://xiucheng.org/assets/pdfs/nips24-sumba.pdf) | [Link](https://github.com/chenxiaodanhit/Sumba/) | NeurIPS'24 | LTSF |
| GLAFF | Rethinking the Power of Timestamps for Robust Time Series Forecasting: A Global-Local Fusion Perspective | [Link](https://arxiv.org/pdf/2409.18696) | [Link](https://github.com/ForestsKing/GLAFF) | NeurIPS'24 | LTSF |
Expand Down
1 change: 1 addition & 0 deletions README_CN.md
Original file line number Diff line number Diff line change
Expand Up @@ -135,6 +135,7 @@ BasicTS 实现了丰富的基线模型,包括经典模型、时空预测模型
| 📊Baseline | 📝Title | 📄Paper | 💻Code | 🏛Venue | 🎯Task |
| :------------ | :------------------------------------------------------------------------------------------------------- | :----------------------------------------------------- | :---------------------------------------------------------------------------- | :--------- | :----- |
| SOFTS | SOFTS: Efficient Multivariate Time Series Forecasting with Series-Core Fusion | [Link](https://arxiv.org/pdf/2404.14197) | [Link](https://github.com/Secilia-Cxy/SOFTS) | NeurIPS'24 | LTSF |
| CATS | Are Self-Attentions Effective for Time Series Forecasting? | [Link](https://arxiv.org/pdf/2405.16877) | [Link](https://github.com/dongbeank/CATS) | NeurIPS'24 | LTSF |
| Sumba | Structured Matrix Basis for Multivariate Time Series Forecasting with Interpretable Dynamics | [Link](https://xiucheng.org/assets/pdfs/nips24-sumba.pdf) | [Link](https://github.com/chenxiaodanhit/Sumba/) | NeurIPS'24 | LTSF |
| GLAFF | Rethinking the Power of Timestamps for Robust Time Series Forecasting: A Global-Local Fusion Perspective | [Link](https://arxiv.org/pdf/2409.18696) | [Link](https://github.com/ForestsKing/GLAFF) | NeurIPS'24 | LTSF |
Expand Down
149 changes: 149 additions & 0 deletions baselines/SOFTS/ETTh1.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,149 @@
import os
import sys
from easydict import EasyDict
sys.path.append(os.path.abspath(__file__ + '/../../..'))
from basicts.metrics import masked_mae, masked_mse, masked_mape, masked_rmse
from basicts.data import TimeSeriesForecastingDataset
from basicts.runners import SimpleTimeSeriesForecastingRunner
from basicts.scaler import ZScoreScaler
from basicts.utils import get_regular_settings

from .arch import SOFTS

############################## Hot Parameters ##############################
# Dataset & Metrics configuration
DATA_NAME = 'ETTh1' # Dataset name
regular_settings = get_regular_settings(DATA_NAME)
INPUT_LEN = regular_settings['INPUT_LEN'] # 336, better performance
OUTPUT_LEN = regular_settings['OUTPUT_LEN'] # Length of output sequence
TRAIN_VAL_TEST_RATIO = regular_settings['TRAIN_VAL_TEST_RATIO'] # Train/Validation/Test split ratios
NORM_EACH_CHANNEL = regular_settings['NORM_EACH_CHANNEL'] # Whether to normalize each channel of the data
RESCALE = regular_settings['RESCALE'] # Whether to rescale the data
NULL_VAL = regular_settings['NULL_VAL'] # Null value in the data
# Model architecture and parameters
MODEL_ARCH = SOFTS
NUM_NODES = 7
MODEL_PARAM = {
"enc_in": NUM_NODES, # num nodes
"dec_in": NUM_NODES,
"c_out": NUM_NODES,
"seq_len": INPUT_LEN,
"pred_len": OUTPUT_LEN, # prediction sequence length
"e_layers": 2, # num of encoder layers
"d_model": 256,
"d_core": 256,
"d_ff": 512,
"dropout": 0.0,
"use_norm" : True,
"activation": "gelu",
"num_time_features": 4, # number of used time features
"time_of_day_size": 24,
"day_of_week_size": 7,
"day_of_month_size": 31,
"day_of_year_size": 366
}
NUM_EPOCHS = 50

############################## General Configuration ##############################
CFG = EasyDict()
# General settings
CFG.DESCRIPTION = 'An Example Config'
CFG.GPU_NUM = 1 # Number of GPUs to use (0 for CPU mode)
# Runner
CFG.RUNNER = SimpleTimeSeriesForecastingRunner

############################## Dataset Configuration ##############################
CFG.DATASET = EasyDict()
# Dataset settings
CFG.DATASET.NAME = DATA_NAME
CFG.DATASET.TYPE = TimeSeriesForecastingDataset
CFG.DATASET.PARAM = EasyDict({
'dataset_name': DATA_NAME,
'train_val_test_ratio': TRAIN_VAL_TEST_RATIO,
'input_len': INPUT_LEN,
'output_len': OUTPUT_LEN,
# 'mode' is automatically set by the runner
})

############################## Scaler Configuration ##############################
CFG.SCALER = EasyDict()
# Scaler settings
CFG.SCALER.TYPE = ZScoreScaler # Scaler class
CFG.SCALER.PARAM = EasyDict({
'dataset_name': DATA_NAME,
'train_ratio': TRAIN_VAL_TEST_RATIO[0],
'norm_each_channel': NORM_EACH_CHANNEL,
'rescale': RESCALE,
})

############################## Model Configuration ##############################
CFG.MODEL = EasyDict()
# Model settings
CFG.MODEL.NAME = MODEL_ARCH.__name__
CFG.MODEL.ARCH = MODEL_ARCH
CFG.MODEL.PARAM = MODEL_PARAM
CFG.MODEL.FORWARD_FEATURES = [0, 1, 2, 3, 4]
CFG.MODEL.TARGET_FEATURES = [0]

############################## Metrics Configuration ##############################

CFG.METRICS = EasyDict()
# Metrics settings
CFG.METRICS.FUNCS = EasyDict({
'MAE': masked_mae,
'MSE': masked_mse,
'RMSE': masked_rmse,
'MAPE': masked_mape
})
CFG.METRICS.TARGET = 'MAE'
CFG.METRICS.NULL_VAL = NULL_VAL

############################## Training Configuration ##############################
CFG.TRAIN = EasyDict()
CFG.TRAIN.NUM_EPOCHS = NUM_EPOCHS
CFG.TRAIN.CKPT_SAVE_DIR = os.path.join(
'checkpoints',
MODEL_ARCH.__name__,
'_'.join([DATA_NAME, str(CFG.TRAIN.NUM_EPOCHS), str(INPUT_LEN), str(OUTPUT_LEN)])
)
CFG.TRAIN.LOSS = masked_mae
# Optimizer settings
CFG.TRAIN.OPTIM = EasyDict()
CFG.TRAIN.OPTIM.TYPE = "Adam"
CFG.TRAIN.OPTIM.PARAM = {
"lr": 0.0003,
}
# Learning rate scheduler settings
CFG.TRAIN.LR_SCHEDULER = EasyDict()
CFG.TRAIN.LR_SCHEDULER.TYPE = "MultiStepLR"
CFG.TRAIN.LR_SCHEDULER.PARAM = {
"milestones": [1, 25, 50],
"gamma": 0.5
}
CFG.TRAIN.CLIP_GRAD_PARAM = {
'max_norm': 5.0
}
# Train data loader settings
CFG.TRAIN.DATA = EasyDict()
CFG.TRAIN.DATA.BATCH_SIZE = 64
CFG.TRAIN.DATA.SHUFFLE = True
CFG.TRAIN.EARLY_STOPPING_PATIENCE = 10

############################## Validation Configuration ##############################
CFG.VAL = EasyDict()
CFG.VAL.INTERVAL = 1
CFG.VAL.DATA = EasyDict()
CFG.VAL.DATA.BATCH_SIZE = 64

############################## Test Configuration ##############################
CFG.TEST = EasyDict()
CFG.TEST.INTERVAL = 1
CFG.TEST.DATA = EasyDict()
CFG.TEST.DATA.BATCH_SIZE = 64

############################## Evaluation Configuration ##############################

CFG.EVAL = EasyDict()

# Evaluation parameters
CFG.EVAL.USE_GPU = True # Whether to use GPU for evaluation. Default: True
153 changes: 153 additions & 0 deletions baselines/SOFTS/ETTh2.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,153 @@
import os
import sys
from easydict import EasyDict
sys.path.append(os.path.abspath(__file__ + '/../../..'))
from basicts.metrics import masked_mae, masked_mse, masked_mape, masked_rmse
from basicts.data import TimeSeriesForecastingDataset
from basicts.runners import SimpleTimeSeriesForecastingRunner
from basicts.scaler import ZScoreScaler
from basicts.utils import get_regular_settings

from .arch import SOFTS

############################## Hot Parameters ##############################
# Dataset & Metrics configuration
DATA_NAME = 'ETTh2' # Dataset name
regular_settings = get_regular_settings(DATA_NAME)
# INPUT_LEN = regular_settings['INPUT_LEN'] # Length of input sequence
INPUT_LEN = 192 # better performance
OUTPUT_LEN = regular_settings['OUTPUT_LEN'] # Length of output sequence
TRAIN_VAL_TEST_RATIO = regular_settings['TRAIN_VAL_TEST_RATIO'] # Train/Validation/Test split ratios
NORM_EACH_CHANNEL = regular_settings['NORM_EACH_CHANNEL'] # Whether to normalize each channel of the data
RESCALE = regular_settings['RESCALE'] # Whether to rescale the data
NULL_VAL = regular_settings['NULL_VAL'] # Null value in the data
# Model architecture and parameters
MODEL_ARCH = SOFTS
NUM_NODES = 7
MODEL_PARAM = {
"enc_in": NUM_NODES, # num nodes
"dec_in": NUM_NODES,
"c_out": NUM_NODES,
"seq_len": INPUT_LEN,
"pred_len": OUTPUT_LEN, # prediction sequence length
"e_layers": 2, # num of encoder layers
"d_model": 128,
"d_core": 64,
"d_ff": 128,
"dropout": 0.0,
"use_norm" : True,
"activation": "gelu",
"num_time_features": 4, # number of used time features
"time_of_day_size": 24,
"day_of_week_size": 7,
"day_of_month_size": 31,
"day_of_year_size": 366
}
NUM_EPOCHS = 20

############################## General Configuration ##############################
CFG = EasyDict()
# General settings
CFG.DESCRIPTION = 'An Example Config'
CFG.GPU_NUM = 1 # Number of GPUs to use (0 for CPU mode)
# Runner
CFG.RUNNER = SimpleTimeSeriesForecastingRunner

CFG.ENV = EasyDict() # Environment settings. Default: None
CFG.ENV.SEED = 2024 # Random seed. Default: None

############################## Dataset Configuration ##############################
CFG.DATASET = EasyDict()
# Dataset settings
CFG.DATASET.NAME = DATA_NAME
CFG.DATASET.TYPE = TimeSeriesForecastingDataset
CFG.DATASET.PARAM = EasyDict({
'dataset_name': DATA_NAME,
'train_val_test_ratio': TRAIN_VAL_TEST_RATIO,
'input_len': INPUT_LEN,
'output_len': OUTPUT_LEN,
# 'mode' is automatically set by the runner
})

############################## Scaler Configuration ##############################
CFG.SCALER = EasyDict()
# Scaler settings
CFG.SCALER.TYPE = ZScoreScaler # Scaler class
CFG.SCALER.PARAM = EasyDict({
'dataset_name': DATA_NAME,
'train_ratio': TRAIN_VAL_TEST_RATIO[0],
'norm_each_channel': NORM_EACH_CHANNEL,
'rescale': RESCALE,
})

############################## Model Configuration ##############################
CFG.MODEL = EasyDict()
# Model settings
CFG.MODEL.NAME = MODEL_ARCH.__name__
CFG.MODEL.ARCH = MODEL_ARCH
CFG.MODEL.PARAM = MODEL_PARAM
CFG.MODEL.FORWARD_FEATURES = [0, 1, 2, 3, 4]
CFG.MODEL.TARGET_FEATURES = [0]

############################## Metrics Configuration ##############################

CFG.METRICS = EasyDict()
# Metrics settings
CFG.METRICS.FUNCS = EasyDict({
'MAE': masked_mae,
'MSE': masked_mse,
'RMSE': masked_rmse,
'MAPE': masked_mape
})
CFG.METRICS.TARGET = 'MAE'
CFG.METRICS.NULL_VAL = NULL_VAL

############################## Training Configuration ##############################
CFG.TRAIN = EasyDict()
CFG.TRAIN.NUM_EPOCHS = NUM_EPOCHS
CFG.TRAIN.CKPT_SAVE_DIR = os.path.join(
'checkpoints',
MODEL_ARCH.__name__,
'_'.join([DATA_NAME, str(CFG.TRAIN.NUM_EPOCHS), str(INPUT_LEN), str(OUTPUT_LEN)])
)
CFG.TRAIN.LOSS = masked_mse
# Optimizer settings
CFG.TRAIN.OPTIM = EasyDict()
CFG.TRAIN.OPTIM.TYPE = "Adam"
CFG.TRAIN.OPTIM.PARAM = {
"lr": 0.0003,
}
# Learning rate scheduler settings
CFG.TRAIN.LR_SCHEDULER = EasyDict()
CFG.TRAIN.LR_SCHEDULER.TYPE = "MultiStepLR"
CFG.TRAIN.LR_SCHEDULER.PARAM = {
"milestones": [1, 25, 50],
"gamma": 0.5
}
CFG.TRAIN.CLIP_GRAD_PARAM = {
'max_norm': 5.0
}
# Train data loader settings
CFG.TRAIN.DATA = EasyDict()
CFG.TRAIN.DATA.BATCH_SIZE = 32
CFG.TRAIN.DATA.SHUFFLE = True
CFG.TRAIN.EARLY_STOPPING_PATIENCE = 10

############################## Validation Configuration ##############################
CFG.VAL = EasyDict()
CFG.VAL.INTERVAL = 1
CFG.VAL.DATA = EasyDict()
CFG.VAL.DATA.BATCH_SIZE = 64

############################## Test Configuration ##############################
CFG.TEST = EasyDict()
CFG.TEST.INTERVAL = 1
CFG.TEST.DATA = EasyDict()
CFG.TEST.DATA.BATCH_SIZE = 64

############################## Evaluation Configuration ##############################

CFG.EVAL = EasyDict()

# Evaluation parameters
CFG.EVAL.USE_GPU = True # Whether to use GPU for evaluation. Default: True
2 changes: 1 addition & 1 deletion baselines/SOFTS/Weather.py
Original file line number Diff line number Diff line change
Expand Up @@ -14,7 +14,7 @@
# Dataset & Metrics configuration
DATA_NAME = 'Weather' # Dataset name
regular_settings = get_regular_settings(DATA_NAME)
INPUT_LEN = regular_settings['INPUT_LEN'] # Length of input sequence
INPUT_LEN = regular_settings['INPUT_LEN'] # 336, better performance
OUTPUT_LEN = regular_settings['OUTPUT_LEN'] # Length of output sequence
TRAIN_VAL_TEST_RATIO = regular_settings['TRAIN_VAL_TEST_RATIO'] # Train/Validation/Test split ratios
NORM_EACH_CHANNEL = regular_settings['NORM_EACH_CHANNEL'] # Whether to normalize each channel of the data
Expand Down
11 changes: 2 additions & 9 deletions baselines/SOFTS/arch/softs_arch.py
Original file line number Diff line number Diff line change
Expand Up @@ -54,7 +54,7 @@ class SOFTS(nn.Module):
'''
Paper: SOFTS: Efficient Multivariate Time Series Forecasting with Series-Core Fusion
Official Code: https://github.com/Secilia-Cxy/SOFTS
Link: https://xiucheng.org/assets/pdfs/nips24-sumba.pdf
Link: https://arxiv.org/pdf/2404.14197
Venue: NeurIPS 2024
Task: Long-term Time Series Forecasting
'''
Expand Down Expand Up @@ -119,18 +119,11 @@ def forward(self, history_data: torch.Tensor, future_data: torch.Tensor, batch_s
torch.Tensor: outputs with shape [B, L2, N, 1]
"""

# change MinuteOfDay to MinuteOfHour
# change TimeOfDay to MinuteOfHour
history_data[..., 1] = history_data[..., 1] * self.time_of_day_size // (self.time_of_day_size / 24) / 23.0
x_enc, x_mark_enc, x_dec, x_mark_dec = data_transformation_4_xformer(history_data=history_data,
future_data=future_data,
start_token_len=0)
#print(x_mark_enc.shape, x_mark_dec.shape)
prediction = self.forward_xformer(x_enc=x_enc, x_mark_enc=x_mark_enc)
return prediction.unsqueeze(-1)







0 comments on commit c2940f2

Please sign in to comment.