forked from NVIDIA/Megatron-LM
-
Notifications
You must be signed in to change notification settings - Fork 0
/
pretrain_vlm.py
282 lines (217 loc) · 10.8 KB
/
pretrain_vlm.py
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
# Copyright (c) 2024, NVIDIA CORPORATION. All rights reserved.
"""Pretrain vision language model."""
from copy import deepcopy
from functools import partial
from types import SimpleNamespace
import torch
from megatron.core import parallel_state, tensor_parallel
from megatron.core.datasets.blended_megatron_dataset_builder import BlendedMegatronDatasetBuilder
from megatron.core.datasets.gpt_dataset import MockGPTLowLevelDataset
from megatron.core.datasets.multimodal_dataset import MockMultimodalDataset, MultimodalDatasetConfig
from megatron.core.enums import ModelType
from megatron.core.models.multimodal.llava_model import LLaVAModel
from megatron.core.models.multimodal.llava_spec import decoder_model_with_transformer_engine_default_spec
from megatron.core.models.vision.vit_layer_specs import get_vit_layer_with_transformer_engine_spec
from megatron.core.transformer.spec_utils import import_module
from megatron.training import get_args, get_timers, get_tokenizer, pretrain, print_rank_0
from megatron.training.arguments import core_transformer_config_from_args
from pretrain_gpt import loss_func
def get_num_image_tokens():
args = get_args()
add_class_token = not args.disable_vision_class_token
num_patches_per_dim_h = args.img_h // args.patch_dim
num_patches_per_dim_w = args.img_w // args.patch_dim
num_patches = num_patches_per_dim_h * num_patches_per_dim_w
num_image_tokens = num_patches + (1 if add_class_token else 0)
return num_image_tokens
def model_provider(
pre_process=True, post_process=True, add_encoder=True, add_decoder=True,
parallel_output=True) -> LLaVAModel:
"""Builds the model.
Note: currently, only LLaVA model is supported. Follow-up changes will make this configurable.
Args:
pre_process (bool): Include the embedding layer in the gpt decoder (used with pipeline parallelism). Defaults to True.
post_process (bool): Include an output layer and a layernorm in the gpt decoder (used with pipeline parallelism). Defaults to True.
add_encoder (bool): Construct the encoder module (used with pipeline parallelism). Defaults to True. When we use pipelining, the encoder
will live on only a subset of the pipeline stages (specifically, only the first stage).
add_decoder (bool): Construct the decoder module (used with pipeline parallelism). Defaults to True. When we use pipelining, the decoder
will live on only a subset of the pipeline stages (specifically, every stage after the first one).
parallel_output (bool): Enable model parallel output.
Returns:
model (megatron.core.models.multimodal.llava_model.LLaVAModel): A multimodal model
"""
args = get_args()
num_image_tokens = get_num_image_tokens()
args.decoder_seq_length = args.seq_length + num_image_tokens
args.seq_length = num_image_tokens
args.max_position_embeddings = max(args.max_position_embeddings, args.decoder_seq_length)
print_rank_0('building a multimodal model ...')
language_transformer_config = core_transformer_config_from_args(get_args())
if args.spec is not None:
language_transformer_layer_spec = import_module(args.spec)
else:
language_transformer_layer_spec = decoder_model_with_transformer_engine_default_spec(
args.num_experts, args.moe_grouped_gemm
)
vision_transformer_layer_spec = get_vit_layer_with_transformer_engine_spec()
# TODO: Make these configurable via input .yaml config.
vision_transformer_config = deepcopy(language_transformer_config)
vision_transformer_config.num_layers = args.encoder_num_layers
if args.pipeline_model_parallel_size > 1:
assert args.encoder_pipeline_model_parallel_size == 1, "ViT can only live on 1 pipeline stage."
vision_transformer_config.pipeline_model_parallel_size = args.encoder_pipeline_model_parallel_size
vision_projection_type = "mlp"
vision_projection_config = deepcopy(language_transformer_config)
vision_projection_modules = deepcopy(language_transformer_layer_spec.submodules.mlp.submodules)
model = LLaVAModel(
language_transformer_config=language_transformer_config,
language_transformer_layer_spec=language_transformer_layer_spec,
language_vocab_size=args.padded_vocab_size,
language_max_sequence_length=args.max_position_embeddings,
vision_transformer_config=vision_transformer_config,
vision_transformer_layer_spec=vision_transformer_layer_spec,
drop_vision_class_token=args.disable_vision_class_token,
vision_projection_config=vision_projection_config,
vision_projection_layer_spec=vision_projection_modules,
vision_projection_type=vision_projection_type,
parallel_output=parallel_output,
language_position_embedding_type=args.position_embedding_type,
language_rotary_percent=args.rotary_percent,
pre_process=pre_process,
post_process=post_process,
add_encoder=add_encoder,
add_decoder=add_decoder,
img_h=args.img_h,
img_w=args.img_w,
patch_dim=args.patch_dim,
)
return model
def train_valid_test_datasets_provider(train_val_test_num_samples):
"""Build the train test and validation datasets.
Args:
train_val_test_num_samples : A list containing the number of samples in train, validation, and test sets.
Returns:
train_ds, val_ds, test_ds (megatron.core.datasets.multimodal_dataset.MockMultimodalDataset): Train, validation, and test datasets, respectively.
"""
args = get_args()
config = MultimodalDatasetConfig(
random_seed=args.seed,
split=args.split,
sequence_length=args.decoder_seq_length-args.seq_length,
tokenizer=get_tokenizer(),
reset_position_ids=args.reset_position_ids,
reset_attention_mask=args.reset_attention_mask,
eod_mask_loss=args.eod_mask_loss,
image_h=args.img_h,
image_w=args.img_w,
preprocess_func=_preprocess_data_for_llava,
)
print_rank_0("> building train, validation, and test datasets for multimodal ...")
train_ds, valid_ds, test_ds = BlendedMegatronDatasetBuilder(
MockMultimodalDataset, train_val_test_num_samples,
lambda: parallel_state.get_tensor_model_parallel_rank() == 0, config
).build()
print_rank_0("> finished creating multimodal datasets ...")
return train_ds, valid_ds, test_ds
def _preprocess_data_for_llava(data):
"""Preprocess data sample to the format expected by a LLaVA model.
Note: This doesn't support all the different modes in the official LLaVA repo yet.
Args:
data (dict): Data sample with keys like 'image', 'tokens', etc.
Returns:
data (dict): Processed data sample suitable for the model.
"""
args = get_args()
# TODO: Move these to multimodal spec (added in a separate code change).
num_image_tokens = get_num_image_tokens()
data["loss_mask"] = torch.cat(
[torch.zeros(num_image_tokens, dtype=torch.float32), data["loss_mask"]]
)
data["labels"] = torch.cat([torch.zeros(num_image_tokens, dtype=torch.int64), data["labels"]])
full_seq_length = len(data["labels"])
attention_mask = torch.tril(torch.ones((1, full_seq_length, full_seq_length)))
attention_mask = attention_mask < 0.5
attention_mask[:, num_image_tokens:, num_image_tokens:] = data["attention_mask"]
data["attention_mask"] = attention_mask
return data
def get_batch(data_iterator):
"""Generate a batch.
Args:
data_iterator: Iterable dataset.
Returns:
sample: A data sample with images, tokens, etc.
"""
# Broadcast data.
if data_iterator is not None:
data = next(data_iterator)
else:
data = None
data_i = tensor_parallel.broadcast_data(["tokens", "position_ids", "labels"], data, torch.int64)
data_f = tensor_parallel.broadcast_data(["image", "loss_mask"], data, torch.float32)
data_b = tensor_parallel.broadcast_data(["attention_mask"], data, torch.bool)
tokens = data_i["tokens"].long()
position_ids = data_i["position_ids"].long()
labels = data_i["labels"].long()
images = data_f["image"].float()
loss_mask = data_f["loss_mask"].float()
attention_mask = data_b["attention_mask"].bool()
return tokens, position_ids, labels, images, loss_mask, attention_mask
def forward_step(data_iterator, model: LLaVAModel):
"""Forward training step.
Args:
data_iterator: Iterable dataset.
model (megatron.core.models.multimodal.llava_model.LLaVAModel): Multimodal model
Returns:
output_tensor (torch.Tensor): Loss of shape [b, s] if labels are provided, otherwise logits of shape [b, s, vocab_size].
loss_func (callable): Loss function with a loss mask specified.
"""
timers = get_timers()
# Get the batch.
timers('batch-generator', log_level=2).start()
tokens, position_ids, labels, images, loss_mask, attention_mask = get_batch(data_iterator)
timers('batch-generator').stop()
output_tensor = model(images, tokens, position_ids, attention_mask, labels=labels)
return output_tensor, partial(loss_func, loss_mask)
def add_vlm_extra_args(parser):
"""Extra arguments."""
group = parser.add_argument_group(title='vision language model specific arguments')
group.add_argument("--disable-vision-class-token", action="store_true", default=False)
return parser
def llava_embedding_ranks(pp_ranks):
"""LLava's embedding ranks consist of the decoder's first and last ranks (ie, the ViT has no embeddings).
Args:
pp_ranks: A list of global ranks that constitute a pipeline group.
"""
args = get_args()
# encoder size is also the index to the first rank of the decoder.
epp = args.encoder_pipeline_model_parallel_size
last_rank = pp_ranks[-1]
if len(pp_ranks) == 1 or pp_ranks[epp] == last_rank:
return [last_rank]
else:
return [pp_ranks[epp], last_rank]
def llava_position_embedding_ranks(pp_ranks):
"""LLava's embedding ranks consist of the singular rank of the model or the decoder's first rank.
Args:
pp_ranks: A list of global ranks that constitute a pipeline group.
"""
args = get_args()
# encoder size is also the index to the first rank of the decoder.
epp = args.encoder_pipeline_model_parallel_size
last_rank = pp_ranks[-1]
if len(pp_ranks) == 1:
return [last_rank]
else:
return [pp_ranks[epp]]
if __name__ == "__main__":
train_valid_test_datasets_provider.is_distributed = True
pretrain(
train_valid_test_datasets_provider,
model_provider,
ModelType.encoder_and_decoder,
forward_step,
args_defaults={'tokenizer_type': 'GPT2BPETokenizer'},
extra_args_provider=add_vlm_extra_args,
get_embedding_ranks=llava_embedding_ranks,
get_position_embedding_ranks=llava_position_embedding_ranks,
)