This is an annotated PyTorch experiment to train a switch transformer.
14import torch
15import torch.nn as nn
16
17from labml import experiment, tracker
18from labml.configs import option
19from labml_helpers.module import Module
20from labml_helpers.train_valid import BatchIndex
21from labml_nn.experiments.nlp_autoregression import NLPAutoRegressionConfigs
24class AutoregressiveModel(Module):
29 def __init__(self, n_vocab: int, d_model: int, transformer: Module):
30 super().__init__()
Token embedding module
32 self.src_embed = nn.Embedding(n_vocab, d_model)
Transformer
34 self.transformer = transformer
Final layer
36 self.generator = nn.Linear(d_model, n_vocab)
37 self.mask = None
39 def forward(self, x: torch.Tensor):
Initialize the subsequent mask
41 if self.mask is None or self.mask.size(0) != len(x):
42 from labml_nn.transformers.utils import subsequent_mask
43 self.mask = subsequent_mask(len(x)).to(x.device)
Token embeddings
45 x = self.src_embed(x)
Run it through the transformer
47 res, counts, route_prob, n_dropped, route_prob_max = self.transformer(x, self.mask)
Generate logits of the next token
49 res = self.generator(res)
51 return res, counts, route_prob, n_dropped, route_prob_max
This extends NLPAutoRegressionConfigs
.
The default configs can and will be over-ridden when we start the experiment
54class Configs(NLPAutoRegressionConfigs):
63 model: AutoregressiveModel
64 transformer: Module
Token embedding size
67 d_model: int = 128
Number of attention heads
69 heads: int = 4
Dropout probability
71 dropout: float = 0.0
Number of features in FFN hidden layer
73 d_ff: int = 256
Number of transformer layers
75 n_layers: int = 6
Number of experts
77 n_experts: int = 4
Load balancing coefficient
79 load_balancing_loss_ceof = 0.01
Whether to scale the chosen expert outputs by the routing probability
81 is_scale_prob: bool = True
Whether to drop tokens
83 drop_tokens: bool = False
Capacity factor to determine capacity of each model
85 capacity_factor: float = 1.0
87 def init(self):
88 super().init()
Initialize tracking indicators
90 tracker.set_scalar("lb_loss.*", False)
91 tracker.set_scalar("route.*", False)
92 tracker.set_scalar("dropped.*", False)
94 def step(self, batch: any, batch_idx: BatchIndex):
Move data to the device
100 data, target = batch[0].to(self.device), batch[1].to(self.device)
Update global step (number of tokens processed) when in training mode
103 if self.mode.is_train:
104 tracker.add_global_step(data.shape[0] * data.shape[1])
Whether to capture model outputs
107 with self.mode.update(is_log_activations=batch_idx.is_last):
Get model outputs.
109 output, counts, route_prob, n_dropped, route_prob_max = self.model(data)
Calculate and cross entropy loss
112 cross_entropy_loss = self.loss_func(output, target)
Total number of tokens processed, , in the current batch
114 total = counts.sum(dim=-1, keepdims=True)
Fraction of tokens routed to each expert is the count of tokens where the argmax of is equal to .
118 route_frac = counts / total
Mean routing probability
121 route_prob = route_prob / total
Load balancing loss is the loss for a single layer and here we are taking the sum of losses across all layers.
126 load_balancing_loss = self.n_experts * (route_frac * route_prob).sum()
Track stats
129 tracker.add('dropped.', total.new_tensor(n_dropped) / total)
130 tracker.add('route.min.', route_frac.min())
131 tracker.add('route.max.', route_frac.max())
132 tracker.add('route.std.', route_frac.std())
133 tracker.add('route.max_prob.', route_prob_max)
134 tracker.add("loss.", cross_entropy_loss)
135 tracker.add("lb_loss.", load_balancing_loss)
Combined loss. The load balancing loss is multiplied by a coefficient which is set to something small like .
140 loss = cross_entropy_loss + self.load_balancing_loss_ceof * load_balancing_loss
Calculate and log accuracy
143 self.accuracy(output, target)
144 self.accuracy.track()
Train the model
147 if self.mode.is_train:
Calculate gradients
149 loss.backward()
Clip gradients
151 torch.nn.utils.clip_grad_norm_(self.model.parameters(), max_norm=self.grad_norm_clip)
Take optimizer step
153 self.optimizer.step()
Log the model parameters and gradients on last batch of every epoch
155 if batch_idx.is_last:
156 tracker.add('model', self.model)
Clear the gradients
158 self.optimizer.zero_grad()
Save the tracked metrics
161 tracker.save()
164@option(Configs.model)
165def autoregressive_model(c: Configs):
169 m = AutoregressiveModel(c.n_tokens, c.d_model, c.transformer)
170 return m.to(c.device)
173@option(Configs.transformer)
174def switch_transformer(c: Configs):
178 from labml_nn.transformers.switch import SwitchTransformer, SwitchTransformerLayer, SwitchFeedForward
179 from labml_nn.transformers import MultiHeadAttention
180 from labml_nn.transformers.feed_forward import FeedForward
181
182 return SwitchTransformer(
183 SwitchTransformerLayer(d_model=c.d_model,
184 attn=MultiHeadAttention(c.heads, c.d_model, c.dropout),
185 feed_forward=SwitchFeedForward(capacity_factor=c.capacity_factor,
186 drop_tokens=c.drop_tokens,
187 is_scale_prob=c.is_scale_prob,
188 n_experts=c.n_experts,
189 expert=FeedForward(c.d_model, c.d_ff, c.dropout),
190 d_model=c.d_model),
191 dropout_prob=c.dropout),
192 c.n_layers)
195def main():
Create experiment
200 experiment.create(name="switch_transformer", comment='')
Create configs
202 conf = Configs()
Load configurations
204 experiment.configs(conf,
A dictionary of configurations to override
206 {'tokenizer': 'character',
207 'text': 'tiny_shakespeare',
208 'optimizer.learning_rate': 1.,
209 'optimizer.optimizer': 'Noam',
210 'prompt': 'It is',
211 'prompt_separator': '',
212
213 'transformer': 'switch_transformer',
214 'n_experts': 4,
215
216 'drop_tokens': True,
217 'capacity_factor': 1.2,
218
219 'train_loader': 'shuffled_train_loader',
220 'valid_loader': 'shuffled_valid_loader',
221
222 'seq_len': 64,
223 'epochs': 128,
224 'batch_size': 32,
225 'inner_iterations': 25,
226 })
Set models for saving and loading
229 experiment.add_pytorch_models({'model': conf.model})
Start the experiment
232 with experiment.start():
TrainValidConfigs.run
234 conf.run()
238if __name__ == '__main__':
239 main()