add pl_glue example test (#6034)

* add pl_glue example test

* for now just test that it runs, next validate results of eval or predict?

* complete the run_pl_glue test to validate the actual outcome

* worked on my machine, CI gets less accuracy - trying higher epochs

* match run_pl.sh hparms

* more epochs?

* trying higher lr

* for now just test that the script runs to a completion

* correct the comment

* if cuda is available, add --fp16 --gpus=1 to cover more bases

* style
This commit is contained in:
Stas Bekman 2020-08-11 00:16:52 -07:00 committed by GitHub
parent b25cec13c5
commit f6c0680d36
No known key found for this signature in database
GPG Key ID: 4AEE18F83AFDEB23
2 changed files with 44 additions and 2 deletions

View File

@ -21,6 +21,8 @@ import sys
import unittest
from unittest.mock import patch
import torch
SRC_DIRS = [
os.path.join(os.path.dirname(__file__), dirname)
@ -32,6 +34,7 @@ sys.path.extend(SRC_DIRS)
if SRC_DIRS is not None:
import run_generation
import run_glue
import run_pl_glue
import run_language_modeling
import run_squad
@ -76,6 +79,41 @@ class ExamplesTests(unittest.TestCase):
for value in result.values():
self.assertGreaterEqual(value, 0.75)
def test_run_pl_glue(self):
stream_handler = logging.StreamHandler(sys.stdout)
logger.addHandler(stream_handler)
testargs = """
run_pl_glue.py
--model_name_or_path bert-base-cased
--data_dir ./tests/fixtures/tests_samples/MRPC/
--task mrpc
--do_train
--do_predict
--output_dir ./tests/fixtures/tests_samples/temp_dir
--train_batch_size=32
--learning_rate=1e-4
--num_train_epochs=1
--seed=42
--max_seq_length=128
""".split()
if torch.cuda.is_available():
testargs += ["--fp16", "--gpus=1"]
with patch.object(sys, "argv", testargs):
result = run_pl_glue.main()
# for now just testing that the script can run to a completion
self.assertGreater(result["acc"], 0.25)
#
# TODO: this fails on CI - doesn't get acc/f1>=0.75:
#
# # remove all the various *loss* attributes
# result = {k: v for k, v in result.items() if "loss" not in k}
# for k, v in result.items():
# self.assertGreaterEqual(v, 0.75, f"({k})")
#
def test_run_language_modeling(self):
stream_handler = logging.StreamHandler(sys.stdout)
logger.addHandler(stream_handler)

View File

@ -176,7 +176,7 @@ class GLUETransformer(BaseTransformer):
return parser
if __name__ == "__main__":
def main():
parser = argparse.ArgumentParser()
add_generic_args(parser, os.getcwd())
parser = GLUETransformer.add_model_specific_args(parser, os.getcwd())
@ -194,4 +194,8 @@ if __name__ == "__main__":
if args.do_predict:
checkpoints = list(sorted(glob.glob(os.path.join(args.output_dir, "checkpointepoch=*.ckpt"), recursive=True)))
model = model.load_from_checkpoint(checkpoints[-1])
trainer.test(model)
return trainer.test(model)
if __name__ == "__main__":
main()