test_cinn_prim_gelu.py 3.9 KB
Newer Older
G
GGBond8488 已提交
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
# Copyright (c) 2023 PaddlePaddle Authors. All Rights Reserved.
#
# 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.

import unittest

import numpy as np

import paddle
import paddle.nn.functional as F
from paddle.fluid import core

TOLERANCE = {
    "float16": {"rtol": 1e-3, "atol": 1e-3},
    "float32": {"rtol": 1e-6, "atol": 1e-6},
    "float64": {"rtol": 1e-15, "atol": 1e-15},
}

approximate_conds = [True, False]


def apply_to_static(net, use_cinn):
    build_strategy = paddle.static.BuildStrategy()
    build_strategy.build_cinn_pass = use_cinn
    return paddle.jit.to_static(net, build_strategy=build_strategy)


def generate_data(shape, dtype="float32"):
    np_data = np.random.random(shape).astype(dtype)
    return np_data


class PrimeNet(paddle.nn.Layer):
    def __init__(self, approximate):
45
        super().__init__()
G
GGBond8488 已提交
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
        self.fc = paddle.nn.Linear(4, 4)
        self.approximate = approximate

    def forward(self, x):
        # y = self.fc(x)
        out = F.gelu(x, approximate=self.approximate)
        return out


class TestPrimForwardAndBackward(unittest.TestCase):
    """
    Test PrimeNet with @to_static + prim forward + prim backward + cinn v.s Dygraph
    """

    def setUp(self):
        paddle.seed(2022)
        self.shapes = [[2, 4], [64, 16, 4]]
        self.dtypes = ["float16", "float32"]

    def train(self, use_prim, data):
        for approximate in approximate_conds:
            return self._train(use_prim, approximate, data)

    def _train(self, use_prim, approximate, data):
        paddle.seed(2022)
        net = PrimeNet(approximate)
        sgd = paddle.optimizer.SGD(
            learning_rate=0.1, parameters=net.parameters()
        )
        core._set_prim_all_enabled(use_prim)
        if use_prim:
            net = apply_to_static(net, use_prim)

        res = []
X
xiongkun 已提交
80
        self.x = data
G
GGBond8488 已提交
81 82 83 84 85 86 87 88 89 90 91 92 93 94 95
        for _ in range(10):
            out = net(data)
            loss = paddle.mean(out)
            loss.backward()
            sgd.step()
            sgd.clear_grad()

            res.append(out.numpy())
        self.check_prim(net, use_prim)

        return res

    def check_prim(self, net, use_prim):
        if not use_prim:
            return
96 97
        # Please use PartialProgramLayer(second output parameter of get_concrete_program) rather than
        # main_program here, as main_program is original program before to_prim.
X
xiongkun 已提交
98 99 100 101 102 103
        fwd_ops = [
            op.type
            for op in net.forward.get_concrete_program(self.x)[1]
            .train_program.block(0)
            .ops
        ]
G
GGBond8488 已提交
104 105 106 107
        # Ensure that gelu is splitted into small ops
        self.assertTrue('gelu' not in fwd_ops)

    def test_cinn_prim(self):
108 109 110 111 112 113 114 115 116 117 118 119 120 121 122 123 124
        for shape in self.shapes:
            for dtype in self.dtypes:
                if paddle.device.get_device() == "cpu" and dtype == "float16":
                    print("need pass this case")
                    continue
                data = generate_data(shape, dtype)
                data_t = paddle.to_tensor(data)
                data_t.stop_gradient = False
                dy_res = self.train(use_prim=False, data=data_t)
                cinn_res = self.train(use_prim=True, data=data_t)
                for i in range(len(dy_res)):
                    np.testing.assert_allclose(
                        cinn_res[i],
                        dy_res[i],
                        rtol=TOLERANCE[dtype]['rtol'],
                        atol=TOLERANCE[dtype]['atol'],
                    )
G
GGBond8488 已提交
125 126 127 128


if __name__ == '__main__':
    unittest.main()