finetune.py 16.9 KB
Newer Older
1
2
3
4
5
import argparse
import glob
import logging
import os
import time
6
import warnings
7
from collections import defaultdict
8
9
from pathlib import Path
from typing import Dict, List, Tuple
10

11
12
import numpy as np
import pytorch_lightning as pl
13
import torch
14
from packaging import version
15
16
from torch.utils.data import DataLoader

17
from lightning_base import BaseTransformer, add_generic_args, generic_train
18
19
from transformers import MBartTokenizer, T5ForConditionalGeneration
from transformers.modeling_bart import shift_tokens_right
20
21
22


try:
23
    from .callbacks import Seq2SeqLoggingCallback, get_checkpoint_callback, get_early_stopping_callback
24
    from .utils import (
25
        ROUGE_KEYS,
26
        LegacySeq2SeqDataset,
27
        Seq2SeqDataset,
28
        assert_all_frozen,
29
        calculate_bleu,
30
        calculate_rouge,
31
32
33
        flatten_list,
        freeze_params,
        get_git_info,
34
        label_smoothed_nll_loss,
35
36
37
38
39
        lmap,
        pickle_save,
        save_git_info,
        save_json,
        use_task_specific_params,
40
    )
41
except ImportError:
42
    from callbacks import Seq2SeqLoggingCallback, get_checkpoint_callback, get_early_stopping_callback
43
    from utils import (
44
        ROUGE_KEYS,
45
        LegacySeq2SeqDataset,
46
47
        Seq2SeqDataset,
        assert_all_frozen,
48
        calculate_bleu,
49
        calculate_rouge,
50
51
52
        flatten_list,
        freeze_params,
        get_git_info,
53
        label_smoothed_nll_loss,
54
55
56
57
58
        lmap,
        pickle_save,
        save_git_info,
        save_json,
        use_task_specific_params,
59
    )
60
61
62
63

logger = logging.getLogger(__name__)


64
65
66
class SummarizationModule(BaseTransformer):
    mode = "summarization"
    loss_names = ["loss"]
67
    metric_names = ROUGE_KEYS
68
    default_val_metric = "rouge2"
69

70
71
72
73
    def __init__(self, hparams, **kwargs):
        super().__init__(hparams, num_labels=None, mode=self.mode, **kwargs)
        use_task_specific_params(self.model, "summarization")
        save_git_info(self.hparams.output_dir)
74
        self.metrics_save_path = Path(self.output_dir) / "metrics.json"
75
        self.hparams_save_path = Path(self.output_dir) / "hparams.pkl"
76
        pickle_save(self.hparams, self.hparams_save_path)
77
        self.step_count = 0
78
        self.metrics = defaultdict(list)
79

80
81
82
        self.dataset_kwargs: dict = dict(
            data_dir=self.hparams.data_dir,
            max_source_length=self.hparams.max_source_length,
83
            prefix=self.model.config.prefix or "",
84
        )
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
        n_observations_per_split = {
            "train": self.hparams.n_train,
            "val": self.hparams.n_val,
            "test": self.hparams.n_test,
        }
        self.n_obs = {k: v if v >= 0 else None for k, v in n_observations_per_split.items()}

        self.target_lens = {
            "train": self.hparams.max_target_length,
            "val": self.hparams.val_max_target_length,
            "test": self.hparams.test_max_target_length,
        }
        assert self.target_lens["train"] <= self.target_lens["val"], f"target_lens: {self.target_lens}"
        assert self.target_lens["train"] <= self.target_lens["test"], f"target_lens: {self.target_lens}"

        if self.hparams.freeze_embeds:
            self.freeze_embeds()
        if self.hparams.freeze_encoder:
103
104
105
            freeze_params(self.model.get_encoder())
            assert_all_frozen(self.model.get_encoder())

106
        self.hparams.git_sha = get_git_info()["repo_sha"]
107
        self.num_workers = hparams.num_workers
108
        self.decoder_start_token_id = None  # default to config
109
110
111
        if self.model.config.decoder_start_token_id is None and isinstance(self.tokenizer, MBartTokenizer):
            self.decoder_start_token_id = self.tokenizer.lang_code_to_id[hparams.tgt_lang]
            self.model.config.decoder_start_token_id = self.decoder_start_token_id
112
113
114
        self.dataset_class = (
            Seq2SeqDataset if hasattr(self.tokenizer, "prepare_seq2seq_batch") else LegacySeq2SeqDataset
        )
115
116
117
        self.eval_beams = self.model.config.num_beams if self.hparams.eval_beams is None else self.hparams.eval_beams
        assert self.eval_beams >= 1, f"got self.eval_beams={self.eval_beams}. Need an integer > 1"
        self.val_metric = self.default_val_metric if self.hparams.val_metric is None else self.hparams.val_metric
118
119
120

    def freeze_embeds(self):
        """Freeze token embeddings and positional embeddings for bart, just token embeddings for t5."""
121
        try:
122
123
124
125
            freeze_params(self.model.model.shared)
            for d in [self.model.model.encoder, self.model.model.decoder]:
                freeze_params(d.embed_positions)
                freeze_params(d.embed_tokens)
126
        except AttributeError:
127
128
129
130
131
132
133
134
135
136
            freeze_params(self.model.shared)
            for d in [self.model.encoder, self.model.decoder]:
                freeze_params(d.embed_tokens)

    def forward(self, input_ids, **kwargs):
        return self.model(input_ids, **kwargs)

    def ids_to_clean_text(self, generated_ids: List[int]):
        gen_text = self.tokenizer.batch_decode(
            generated_ids, skip_special_tokens=True, clean_up_tokenization_spaces=True
137
        )
138
        return lmap(str.strip, gen_text)
139

140
    def _step(self, batch: dict) -> Tuple:
141
        pad_token_id = self.tokenizer.pad_token_id
142
143
        src_ids, src_mask = batch["input_ids"], batch["attention_mask"]
        tgt_ids = batch["labels"]
144
        if isinstance(self.model, T5ForConditionalGeneration):
145
            decoder_input_ids = self.model._shift_right(tgt_ids)
146
        else:
147
            decoder_input_ids = shift_tokens_right(tgt_ids, pad_token_id)
148

149
150
        outputs = self(src_ids, attention_mask=src_mask, decoder_input_ids=decoder_input_ids, use_cache=False)
        lm_logits = outputs[0]
151
        if self.hparams.label_smoothing == 0:
152
            # Same behavior as modeling_bart.py, besides ignoring pad_token_id
153
            ce_loss_fct = torch.nn.CrossEntropyLoss(ignore_index=pad_token_id)
154

155
            assert lm_logits.shape[-1] == self.model.config.vocab_size
156
            loss = ce_loss_fct(lm_logits.view(-1, lm_logits.shape[-1]), tgt_ids.view(-1))
157
        else:
158
            lprobs = torch.nn.functional.log_softmax(lm_logits, dim=-1)
159
            loss, nll_loss = label_smoothed_nll_loss(
160
                lprobs, tgt_ids, self.hparams.label_smoothing, ignore_index=pad_token_id
161
            )
162
163
        return (loss,)

164
165
166
167
    @property
    def pad(self) -> int:
        return self.tokenizer.pad_token_id

168
169
    def training_step(self, batch, batch_idx) -> Dict:
        loss_tensors = self._step(batch)
170

171
        logs = {name: loss for name, loss in zip(self.loss_names, loss_tensors)}
172
        # tokens per batch
173
        logs["tpb"] = batch["input_ids"].ne(self.pad).sum() + batch["labels"].ne(self.pad).sum()
174
175
176
177
178
        return {"loss": loss_tensors[0], "log": logs}

    def validation_step(self, batch, batch_idx) -> Dict:
        return self._generative_step(batch)

179
    def validation_epoch_end(self, outputs, prefix="val") -> Dict:
180
181
182
        self.step_count += 1
        losses = {k: torch.stack([x[k] for x in outputs]).mean() for k in self.loss_names}
        loss = losses["loss"]
183
184
185
186
187
188
189
190
191
192
193
194
        generative_metrics = {
            k: np.array([x[k] for x in outputs]).mean() for k in self.metric_names + ["gen_time", "gen_len"]
        }
        metric_val = (
            generative_metrics[self.val_metric] if self.val_metric in generative_metrics else losses[self.val_metric]
        )
        metric_tensor: torch.FloatTensor = torch.tensor(metric_val).type_as(loss)
        generative_metrics.update({k: v.item() for k, v in losses.items()})
        losses.update(generative_metrics)
        all_metrics = {f"{prefix}_avg_{k}": x for k, x in losses.items()}
        all_metrics["step_count"] = self.step_count
        self.save_metrics(all_metrics, prefix)  # writes to self.metrics_save_path
195
        preds = flatten_list([x["preds"] for x in outputs])
196
197
198
199
200
201
        return {
            "log": all_metrics,
            "preds": preds,
            f"{prefix}_loss": loss,
            f"{prefix}_{self.val_metric}": metric_tensor,
        }
202
203
204
205

    def save_metrics(self, latest_metrics, type_path) -> None:
        self.metrics[type_path].append(latest_metrics)
        save_json(self.metrics, self.metrics_save_path)
206

207
208
    def calc_generative_metrics(self, preds, target) -> Dict:
        return calculate_rouge(preds, target)
209

210
    def _generative_step(self, batch: dict) -> dict:
211
        t0 = time.time()
212
        generated_ids = self.model.generate(
213
214
            batch["input_ids"],
            attention_mask=batch["attention_mask"],
215
216
            use_cache=True,
            decoder_start_token_id=self.decoder_start_token_id,
217
            num_beams=self.eval_beams,
218
        )
219
220
        gen_time = (time.time() - t0) / batch["input_ids"].shape[0]
        preds: List[str] = self.ids_to_clean_text(generated_ids)
221
        target: List[str] = self.ids_to_clean_text(batch["labels"])
222
223
        loss_tensors = self._step(batch)
        base_metrics = {name: loss for name, loss in zip(self.loss_names, loss_tensors)}
224
        rouge: Dict = self.calc_generative_metrics(preds, target)
225
        summ_len = np.mean(lmap(len, generated_ids))
226
        base_metrics.update(gen_time=gen_time, gen_len=summ_len, preds=preds, target=target, **rouge)
227
        return base_metrics
228

229
230
    def test_step(self, batch, batch_idx):
        return self._generative_step(batch)
231
232

    def test_epoch_end(self, outputs):
233
        return self.validation_epoch_end(outputs, prefix="test")
234

235
    def get_dataset(self, type_path) -> Seq2SeqDataset:
236
237
        n_obs = self.n_obs[type_path]
        max_target_length = self.target_lens[type_path]
238
        dataset = self.dataset_class(
239
240
241
242
243
244
245
246
            self.tokenizer,
            type_path=type_path,
            n_obs=n_obs,
            max_target_length=max_target_length,
            **self.dataset_kwargs,
        )
        return dataset

247
    def get_dataloader(self, type_path: str, batch_size: int, shuffle: bool = False) -> DataLoader:
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
        dataset = self.get_dataset(type_path)
        sampler = None
        if self.hparams.sortish_sampler and type_path == "train":
            assert self.hparams.gpus <= 1  # TODO: assert earlier
            sampler = dataset.make_sortish_sampler(batch_size)
            shuffle = False

        dataloader = DataLoader(
            dataset,
            batch_size=batch_size,
            collate_fn=dataset.collate_fn,
            shuffle=shuffle,
            num_workers=self.num_workers,
            sampler=sampler,
        )
263
264
265
        return dataloader

    def train_dataloader(self) -> DataLoader:
266
        dataloader = self.get_dataloader("train", batch_size=self.hparams.train_batch_size, shuffle=True)
267
268
        return dataloader

269
270
    def val_dataloader(self) -> DataLoader:
        return self.get_dataloader("val", batch_size=self.hparams.eval_batch_size)
271

272
273
    def test_dataloader(self) -> DataLoader:
        return self.get_dataloader("test", batch_size=self.hparams.eval_batch_size)
274
275
276
277

    @staticmethod
    def add_model_specific_args(parser, root_dir):
        BaseTransformer.add_model_specific_args(parser, root_dir)
278
        add_generic_args(parser, root_dir)
279
        parser.add_argument(
280
            "--max_source_length",
281
282
283
284
285
            default=1024,
            type=int,
            help="The maximum total input sequence length after tokenization. Sequences longer "
            "than this will be truncated, sequences shorter will be padded.",
        )
286
287
288
289
290
291
292
        parser.add_argument(
            "--max_target_length",
            default=56,
            type=int,
            help="The maximum total input sequence length after tokenization. Sequences longer "
            "than this will be truncated, sequences shorter will be padded.",
        )
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
        parser.add_argument(
            "--val_max_target_length",
            default=142,  # these defaults are optimized for CNNDM. For xsum, see README.md.
            type=int,
            help="The maximum total input sequence length after tokenization. Sequences longer "
            "than this will be truncated, sequences shorter will be padded.",
        )
        parser.add_argument(
            "--test_max_target_length",
            default=142,
            type=int,
            help="The maximum total input sequence length after tokenization. Sequences longer "
            "than this will be truncated, sequences shorter will be padded.",
        )
        parser.add_argument("--freeze_encoder", action="store_true")
        parser.add_argument("--freeze_embeds", action="store_true")
        parser.add_argument("--sortish_sampler", action="store_true", default=False)
310
        parser.add_argument("--logger_name", type=str, choices=["default", "wandb", "wandb_shared"], default="default")
311
312
313
        parser.add_argument("--n_train", type=int, default=-1, required=False, help="# examples. -1 means use all.")
        parser.add_argument("--n_val", type=int, default=500, required=False, help="# examples. -1 means use all.")
        parser.add_argument("--n_test", type=int, default=-1, required=False, help="# examples. -1 means use all.")
314
315
316
        parser.add_argument(
            "--task", type=str, default="summarization", required=False, help="# examples. -1 means use all."
        )
317
        parser.add_argument("--label_smoothing", type=float, default=0.0, required=False)
318
319
        parser.add_argument("--src_lang", type=str, default="", required=False)
        parser.add_argument("--tgt_lang", type=str, default="", required=False)
320
        parser.add_argument("--eval_beams", type=int, default=None, required=False)
321
322
323
        parser.add_argument(
            "--val_metric", type=str, default=None, required=False, choices=["bleu", "rouge2", "loss", None]
        )
324
        parser.add_argument("--save_top_k", type=int, default=1, required=False, help="How many checkpoints to save")
325
326
327
328
329
330
331
        parser.add_argument(
            "--early_stopping_patience",
            type=int,
            default=-1,
            required=False,
            help="-1 means never early stop. early_stopping_patience is measured in validation checks, not epochs. So val_check_interval will effect it.",
        )
332
333
334
        return parser


335
336
337
338
class TranslationModule(SummarizationModule):
    mode = "translation"
    loss_names = ["loss"]
    metric_names = ["bleu"]
339
    default_val_metric = "bleu"
340

341
342
343
344
345
    def __init__(self, hparams, **kwargs):
        super().__init__(hparams, **kwargs)
        self.dataset_kwargs["src_lang"] = hparams.src_lang
        self.dataset_kwargs["tgt_lang"] = hparams.tgt_lang

346
    def calc_generative_metrics(self, preds, target) -> dict:
347
        return calculate_bleu(preds, target)
348
349


350
351
352
353
354
def main(args, model=None) -> SummarizationModule:
    Path(args.output_dir).mkdir(exist_ok=True)
    if len(os.listdir(args.output_dir)) > 3 and args.do_train:
        raise ValueError("Output directory ({}) already exists and is not empty.".format(args.output_dir))
    if model is None:
355
        if "summarization" in args.task:
356
357
358
            model: SummarizationModule = SummarizationModule(args)
        else:
            model: SummarizationModule = TranslationModule(args)
359
360
    if version.parse(torch.__version__) == version.parse("1.6") and args.fp16:
        warnings.warn("FP16 only seems to work with torch 1.5+apex")
361
    dataset = Path(args.data_dir).name
362
    if (
363
        args.logger_name == "default"
364
365
366
367
368
        or args.fast_dev_run
        or str(args.output_dir).startswith("/tmp")
        or str(args.output_dir).startswith("/var")
    ):
        logger = True  # don't pollute wandb logs unnecessarily
369
    elif args.logger_name == "wandb":
370
371
        from pytorch_lightning.loggers import WandbLogger

372
373
        project = os.environ.get("WANDB_PROJECT", dataset)
        logger = WandbLogger(name=model.output_dir.name, project=project)
374

375
    elif args.logger_name == "wandb_shared":
376
377
        from pytorch_lightning.loggers import WandbLogger

378
        logger = WandbLogger(name=model.output_dir.name, project=f"hf_{dataset}")
379
380
381
382
383

    if args.early_stopping_patience >= 0:
        es_callback = get_early_stopping_callback(model.val_metric, args.early_stopping_patience)
    else:
        es_callback = False
384
385

    lower_is_better = args.val_metric == "loss"
386
387
388
389
    trainer: pl.Trainer = generic_train(
        model,
        args,
        logging_callback=Seq2SeqLoggingCallback(),
390
391
392
        checkpoint_callback=get_checkpoint_callback(
            args.output_dir, model.val_metric, args.save_top_k, lower_is_better
        ),
393
        early_stopping_callback=es_callback,
394
395
        logger=logger,
    )
396
    pickle_save(model.hparams, model.output_dir / "hparams.pkl")
397
398
399
400
401
402
403
404
405
    if not args.do_predict:
        return model

    model.hparams.test_checkpoint = ""
    checkpoints = list(sorted(glob.glob(os.path.join(args.output_dir, "*.ckpt"), recursive=True)))
    if checkpoints:
        model.hparams.test_checkpoint = checkpoints[-1]
        trainer.resume_from_checkpoint = checkpoints[-1]
    trainer.logger.log_hyperparams(model.hparams)
406
407
408

    # test() without a model tests using the best checkpoint automatically
    trainer.test()
409
    return model
410
411
412
413


if __name__ == "__main__":
    parser = argparse.ArgumentParser()
414
    parser = pl.Trainer.add_argparse_args(parser)
415
    parser = SummarizationModule.add_model_specific_args(parser, os.getcwd())
416

417
418
419
    args = parser.parse_args()

    main(args)