File size: 2,399 Bytes
74e8f2f
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
 
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
# Copyright 2024 Big Vision Authors.
#
# Licensed under the Apache License, Version 2.0 (the "License");
# you may not use this file except in compliance with the License.
# You may obtain a copy of the License at
#
#     http://www.apache.org/licenses/LICENSE-2.0
#
# Unless required by applicable law or agreed to in writing, software
# distributed under the License is distributed on an "AS IS" BASIS,
# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
# See the License for the specific language governing permissions and
# limitations under the License.

"""Tests for bert."""

import tempfile

from big_vision import input_pipeline
from big_vision.models.proj.flaxformer import bert
from big_vision.models.proj.flaxformer import bert_test_util
import big_vision.pp.builder as pp_builder
import big_vision.pp.ops_general  # pylint: disable=unused-import
import big_vision.pp.proj.flaxformer.bert_ops  # pylint: disable=unused-import
import flax
import jax
import jax.numpy as jnp
import tensorflow as tf


# BERT vocabulary for testing.
_BERT_VOCAB = [
    "[PAD]",
    "[UNK]",
    "this",
    "is",
    "a",
    "test",
    "[CLS]",
    "[SEP]",
]
_TOKEN_LEN = 16


class BertTest(tf.test.TestCase):

  def test_load_apply(self):
    inkey = "text"
    vocab_path = f"{tempfile.mkdtemp()}/vocab.txt"
    with open(vocab_path, "w") as f:
      f.write("\n".join(_BERT_VOCAB))
    ds2, _ = input_pipeline.make_for_inference(
        tf.data.Dataset.from_tensor_slices(
            {inkey: tf.ragged.constant([["this is a test"]])}),
        num_ex_per_process=[1],
        preprocess_fn=pp_builder.get_preprocess_fn(
            f"bert_tokenize(inkey='{inkey}', vocab_path='{vocab_path}', "
            f"max_len={_TOKEN_LEN})"
            "|keep('labels')"),
        batch_size=1,
    )
    text = jnp.array(next(iter(ds2))["labels"])
    model = bert.Model(config="base")
    variables = model.init(jax.random.PRNGKey(0), text)
    params = bert.load(flax.core.unfreeze(variables)["params"],
                       bert_test_util.create_base_checkpoint())
    x, out = model.apply({"params": params}, text)
    self.assertAllEqual(jax.tree_map(jnp.shape, x), (1, 768))
    self.assertAllEqual(
        jax.tree_map(jnp.shape, out), {
            "transformed": (1, 16, 768),
            "pre_logits": (1, 768),
        })


if __name__ == "__main__":
  tf.test.main()