activation.py 8.4 KB
Newer Older
1
"""Custom activation functions."""
2
import math
3
4
from typing import Optional

Woosuk Kwon's avatar
Woosuk Kwon committed
5
6
import torch
import torch.nn as nn
7
import torch.nn.functional as F
Woosuk Kwon's avatar
Woosuk Kwon committed
8

9
10
from vllm.distributed import (divide, get_tensor_model_parallel_rank,
                              get_tensor_model_parallel_world_size)
11
from vllm.model_executor.custom_op import CustomOp
12
from vllm.model_executor.layers.quantization import QuantizationConfig
13
from vllm.model_executor.utils import set_weight_attrs
Woosuk Kwon's avatar
Woosuk Kwon committed
14
15


16
class SiluAndMul(CustomOp):
17
18
    """An activation function for SwiGLU.

19
    The function computes x -> silu(x[:d]) * x[d:] where d = x.shape[-1] // 2.
Woosuk Kwon's avatar
Woosuk Kwon committed
20

21
    Shapes:
22
23
        x: (num_tokens, 2 * d) or (batch_size, seq_len, 2 * d)
        return: (num_tokens, d) or (batch_size, seq_len, d)
24
    """
Woosuk Kwon's avatar
Woosuk Kwon committed
25

26
    def forward_native(self, x: torch.Tensor) -> torch.Tensor:
27
28
29
30
        """PyTorch-native implementation equivalent to forward()."""
        d = x.shape[-1] // 2
        return F.silu(x[..., :d]) * x[..., d:]

31
32
33
    def forward_cuda(self, x: torch.Tensor) -> torch.Tensor:
        from vllm import _custom_ops as ops

34
35
36
        d = x.shape[-1] // 2
        output_shape = (x.shape[:-1] + (d, ))
        out = torch.empty(output_shape, dtype=x.dtype, device=x.device)
37
        ops.silu_and_mul(out, x)
Woosuk Kwon's avatar
Woosuk Kwon committed
38
        return out
39

40
41
42
43
44
45
46
47
48
    def forward_xpu(self, x: torch.Tensor) -> torch.Tensor:
        from vllm._ipex_ops import ipex_ops as ops

        d = x.shape[-1] // 2
        output_shape = (x.shape[:-1] + (d, ))
        out = torch.empty(output_shape, dtype=x.dtype, device=x.device)
        ops.silu_and_mul(out, x)
        return out

49

50
class GeluAndMul(CustomOp):
51
52
53
54
55
56
57
58
59
    """An activation function for GeGLU.

    The function computes x -> GELU(x[:d]) * x[d:] where d = x.shape[-1] // 2.

    Shapes:
        x: (batch_size, seq_len, 2 * d) or (num_tokens, 2 * d)
        return: (batch_size, seq_len, d) or (num_tokens, d)
    """

60
61
62
63
64
65
    def __init__(self, approximate: str = "none"):
        super().__init__()
        self.approximate = approximate
        if approximate not in ("none", "tanh"):
            raise ValueError(f"Unknown approximate mode: {approximate}")

66
    def forward_native(self, x: torch.Tensor) -> torch.Tensor:
67
68
        """PyTorch-native implementation equivalent to forward()."""
        d = x.shape[-1] // 2
69
        return F.gelu(x[..., :d], approximate=self.approximate) * x[..., d:]
70

71
72
73
    def forward_cuda(self, x: torch.Tensor) -> torch.Tensor:
        from vllm import _custom_ops as ops

74
75
76
        d = x.shape[-1] // 2
        output_shape = (x.shape[:-1] + (d, ))
        out = torch.empty(output_shape, dtype=x.dtype, device=x.device)
77
78
79
80
        if self.approximate == "none":
            ops.gelu_and_mul(out, x)
        elif self.approximate == "tanh":
            ops.gelu_tanh_and_mul(out, x)
81
82
        return out

83
84
85
86
87
88
89
90
91
92
93
94
    def forward_xpu(self, x: torch.Tensor) -> torch.Tensor:
        from vllm._ipex_ops import ipex_ops as ops

        d = x.shape[-1] // 2
        output_shape = (x.shape[:-1] + (d, ))
        out = torch.empty(output_shape, dtype=x.dtype, device=x.device)
        if self.approximate == "none":
            ops.gelu_and_mul(out, x)
        elif self.approximate == "tanh":
            ops.gelu_tanh_and_mul(out, x)
        return out

95
96
97
    def extra_repr(self) -> str:
        return f'approximate={repr(self.approximate)}'

98

99
class NewGELU(CustomOp):
100

101
    def forward_native(self, x: torch.Tensor) -> torch.Tensor:
102
103
104
105
106
        """PyTorch-native implementation equivalent to forward()."""
        c = math.sqrt(2.0 / math.pi)
        return 0.5 * x * (1.0 + torch.tanh(c *
                                           (x + 0.044715 * torch.pow(x, 3.0))))

107
108
109
    def forward_cuda(self, x: torch.Tensor) -> torch.Tensor:
        from vllm import _custom_ops as ops

110
        out = torch.empty_like(x)
111
        ops.gelu_new(out, x)
112
113
        return out

114
115
116
    def forward_xpu(self, x: torch.Tensor) -> torch.Tensor:
        from vllm._ipex_ops import ipex_ops as ops

117
        return ops.gelu_new(x)
118

119

120
class FastGELU(CustomOp):
121

122
    def forward_native(self, x: torch.Tensor) -> torch.Tensor:
123
124
125
126
        """PyTorch-native implementation equivalent to forward()."""
        return 0.5 * x * (1.0 + torch.tanh(x * 0.7978845608 *
                                           (1.0 + 0.044715 * x * x)))

127
128
129
    def forward_cuda(self, x: torch.Tensor) -> torch.Tensor:
        from vllm import _custom_ops as ops

130
        out = torch.empty_like(x)
131
        ops.gelu_fast(out, x)
132
133
        return out

134
135
136
    def forward_xpu(self, x: torch.Tensor) -> torch.Tensor:
        from vllm._ipex_ops import ipex_ops as ops

137
        return ops.gelu_fast(x)
138

139

140
141
142
143
144
145
146
147
148
149
150
151
152
153
class QuickGELU(CustomOp):

    # https://github.com/huggingface/transformers/blob/main/src/transformers/activations.py#L90
    def forward_native(self, x: torch.Tensor) -> torch.Tensor:
        """PyTorch-native implementation equivalent to forward()."""
        return x * torch.sigmoid(1.702 * x)

    def forward_cuda(self, x: torch.Tensor) -> torch.Tensor:
        from vllm import _custom_ops as ops

        out = torch.empty_like(x)
        ops.gelu_quick(out, x)
        return out

154
155
156
157
158
159
160
    def forward_xpu(self, x: torch.Tensor) -> torch.Tensor:
        from vllm._ipex_ops import ipex_ops as ops

        out = torch.empty_like(x)
        ops.gelu_quick(out, x)
        return out

161
162
163
    # TODO implement forward_xpu for QuickGELU
    # def forward_xpu(self, x: torch.Tensor) -> torch.Tensor:

164

165
166
167
168
169
170
171
class ReLUSquaredActivation(CustomOp):
    """
    Applies the relu^2 activation introduced in https://arxiv.org/abs/2109.08668v2
    """

    def forward_native(self, x: torch.Tensor) -> torch.Tensor:
        """PyTorch-native implementation equivalent to forward()."""
172
        return torch.square(F.relu(x))
173
174
175
176
177

    def forward_cuda(self, x: torch.Tensor) -> torch.Tensor:
        return self.forward_native(x)


178
179
180
181
182
183
184
185
186
class ScaledActivation(nn.Module):
    """An activation function with post-scale parameters.

    This is used for some quantization methods like AWQ.
    """

    def __init__(
        self,
        act_module: nn.Module,
187
188
189
        intermediate_size: int,
        input_is_parallel: bool = True,
        params_dtype: Optional[torch.dtype] = None,
190
191
192
    ):
        super().__init__()
        self.act = act_module
193
        self.input_is_parallel = input_is_parallel
194
195
196
197
198
199
200
201
        if input_is_parallel:
            tp_size = get_tensor_model_parallel_world_size()
            intermediate_size_per_partition = divide(intermediate_size,
                                                     tp_size)
        else:
            intermediate_size_per_partition = intermediate_size
        if params_dtype is None:
            params_dtype = torch.get_default_dtype()
202
        self.scales = nn.Parameter(
203
            torch.empty(intermediate_size_per_partition, dtype=params_dtype))
204
        set_weight_attrs(self.scales, {"weight_loader": self.weight_loader})
205

206
    def forward(self, x: torch.Tensor) -> torch.Tensor:
207
208
        return self.act(x) / self.scales

209
210
    def weight_loader(self, param: nn.Parameter, loaded_weight: torch.Tensor):
        param_data = param.data
211
212
213
214
215
        if self.input_is_parallel:
            tp_rank = get_tensor_model_parallel_rank()
            shard_size = param_data.shape[0]
            start_idx = tp_rank * shard_size
            loaded_weight = loaded_weight.narrow(0, start_idx, shard_size)
216
217
218
        assert param_data.shape == loaded_weight.shape
        param_data.copy_(loaded_weight)

219

220
221
222
223
224
225
_ACTIVATION_REGISTRY = {
    "gelu": nn.GELU(),
    "gelu_fast": FastGELU(),
    "gelu_new": NewGELU(),
    "gelu_pytorch_tanh": nn.GELU(approximate="tanh"),
    "relu": nn.ReLU(),
226
    "relu2": ReLUSquaredActivation(),
227
    "quick_gelu": QuickGELU(),
228
229
230
}


231
232
233
234
def get_act_fn(
    act_fn_name: str,
    quant_config: Optional[QuantizationConfig] = None,
    intermediate_size: Optional[int] = None,
235
236
    input_is_parallel: bool = True,
    params_dtype: Optional[torch.dtype] = None,
237
) -> nn.Module:
238
    """Get an activation function by name."""
239
240
241
242
243
244
    act_fn_name = act_fn_name.lower()
    if act_fn_name not in _ACTIVATION_REGISTRY:
        raise ValueError(
            f"Activation function {act_fn_name!r} is not supported.")

    act_fn = _ACTIVATION_REGISTRY[act_fn_name]
245
246
    if (quant_config is not None
            and act_fn_name in quant_config.get_scaled_act_names()):
247
248
249
        if intermediate_size is None:
            raise ValueError("intermediate_size must be specified for scaled "
                             "activation functions.")
250
251
        return ScaledActivation(act_fn, intermediate_size, input_is_parallel,
                                params_dtype)
252
    return act_fn