matmul.py 4.54 KB
Newer Older
rusty1s's avatar
rusty1s committed
1
from typing import Tuple
rusty1s's avatar
matmul  
rusty1s committed
2

rusty1s's avatar
rusty1s committed
3
import torch
rusty1s's avatar
rusty1s committed
4

rusty1s's avatar
matmul  
rusty1s committed
5
6
7
8
9
10
11
12
13
14
from torch_sparse.tensor import SparseTensor


def spmm_sum(src: SparseTensor, other: torch.Tensor) -> torch.Tensor:
    rowptr, col, value = src.csr()

    row = src.storage._row
    csr2csc = src.storage._csr2csc
    colptr = src.storage._colptr

rusty1s's avatar
rusty1s committed
15
16
17
    if value is not None:
        value = value.to(other.dtype)

rusty1s's avatar
matmul  
rusty1s committed
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
    if value is not None and value.requires_grad:
        row = src.storage.row()

    if other.requires_grad:
        row = src.storage.row()
        csr2csc = src.storage.csr2csc()
        colptr = src.storage.colptr()

    return torch.ops.torch_sparse.spmm_sum(row, rowptr, col, value, colptr,
                                           csr2csc, other)


def spmm_add(src: SparseTensor, other: torch.Tensor) -> torch.Tensor:
    return spmm_sum(src, other)


34
35
36
37
38
39
40
41
def spmm_mean(src: SparseTensor, other: torch.Tensor) -> torch.Tensor:
    rowptr, col, value = src.csr()

    row = src.storage._row
    rowcount = src.storage._rowcount
    csr2csc = src.storage._csr2csc
    colptr = src.storage._colptr

rusty1s's avatar
rusty1s committed
42
43
44
    if value is not None:
        value = value.to(other.dtype)

45
46
47
48
49
50
51
52
53
54
55
56
57
    if value is not None and value.requires_grad:
        row = src.storage.row()

    if other.requires_grad:
        row = src.storage.row()
        rowcount = src.storage.rowcount()
        csr2csc = src.storage.csr2csc()
        colptr = src.storage.colptr()

    return torch.ops.torch_sparse.spmm_mean(row, rowptr, col, value, rowcount,
                                            colptr, csr2csc, other)


rusty1s's avatar
rusty1s committed
58
59
60
def spmm_min(src: SparseTensor,
             other: torch.Tensor) -> Tuple[torch.Tensor, torch.Tensor]:
    rowptr, col, value = src.csr()
rusty1s's avatar
rusty1s committed
61
62
63
64

    if value is not None:
        value = value.to(other.dtype)

rusty1s's avatar
rusty1s committed
65
66
67
68
69
70
    return torch.ops.torch_sparse.spmm_min(rowptr, col, value, other)


def spmm_max(src: SparseTensor,
             other: torch.Tensor) -> Tuple[torch.Tensor, torch.Tensor]:
    rowptr, col, value = src.csr()
rusty1s's avatar
rusty1s committed
71
72
73
74

    if value is not None:
        value = value.to(other.dtype)

rusty1s's avatar
rusty1s committed
75
76
77
    return torch.ops.torch_sparse.spmm_max(rowptr, col, value, other)


rusty1s's avatar
matmul  
rusty1s committed
78
79
80
81
def spmm(src: SparseTensor, other: torch.Tensor,
         reduce: str = "sum") -> torch.Tensor:
    if reduce == 'sum' or reduce == 'add':
        return spmm_sum(src, other)
82
83
    elif reduce == 'mean':
        return spmm_mean(src, other)
rusty1s's avatar
rusty1s committed
84
85
86
87
88
89
90
91
92
    elif reduce == 'min':
        return spmm_min(src, other)[0]
    elif reduce == 'max':
        return spmm_max(src, other)[0]
    else:
        raise ValueError


def spspmm_sum(src: SparseTensor, other: SparseTensor) -> SparseTensor:
rusty1s's avatar
rusty1s committed
93
    assert src.sparse_size(1) == other.sparse_size(0)
rusty1s's avatar
rusty1s committed
94
95
    rowptrA, colA, valueA = src.csr()
    rowptrB, colB, valueB = other.csr()
rusty1s's avatar
rusty1s committed
96
    value = valueA if valueA is not None else valueB
rusty1s's avatar
rusty1s committed
97
98
    if valueA is not None and valueA.dtype == torch.half:
        valueA = valueA.to(torch.float)
rusty1s's avatar
rusty1s committed
99
100
    if valueB is not None and valueB.dtype == torch.half:
        valueB = valueB.to(torch.float)
rusty1s's avatar
rusty1s committed
101
102
103
    M, K = src.sparse_size(0), other.sparse_size(1)
    rowptrC, colC, valueC = torch.ops.torch_sparse.spspmm_sum(
        rowptrA, colA, valueA, rowptrB, colB, valueB, K)
rusty1s's avatar
rusty1s committed
104
105
    if valueC is not None and value is not None:
        valueC = valueC.to(value.dtype)
106
    return SparseTensor(row=None, rowptr=rowptrC, col=colC, value=valueC,
rusty1s's avatar
rusty1s committed
107
                        sparse_sizes=(M, K), is_sorted=True)
rusty1s's avatar
rusty1s committed
108
109
110
111
112
113
114
115
116
117
118
119


def spspmm_add(src: SparseTensor, other: SparseTensor) -> SparseTensor:
    return spspmm_sum(src, other)


def spspmm(src: SparseTensor, other: SparseTensor,
           reduce: str = "sum") -> SparseTensor:
    if reduce == 'sum' or reduce == 'add':
        return spspmm_sum(src, other)
    elif reduce == 'mean' or reduce == 'min' or reduce == 'max':
        raise NotImplementedError
rusty1s's avatar
matmul  
rusty1s committed
120
121
122
123
    else:
        raise ValueError


rusty1s's avatar
rusty1s committed
124
@torch.jit._overload  # noqa: F811
rusty1s's avatar
rusty1s committed
125
def matmul(src, other, reduce):  # noqa: F811
rusty1s's avatar
rusty1s committed
126
127
128
129
130
    # type: (SparseTensor, torch.Tensor, str) -> torch.Tensor
    pass


@torch.jit._overload  # noqa: F811
rusty1s's avatar
rusty1s committed
131
def matmul(src, other, reduce):  # noqa: F811
rusty1s's avatar
rusty1s committed
132
133
134
135
136
137
    # type: (SparseTensor, SparseTensor, str) -> SparseTensor
    pass


def matmul(src, other, reduce="sum"):  # noqa: F811
    if isinstance(other, torch.Tensor):
rusty1s's avatar
matmul  
rusty1s committed
138
        return spmm(src, other, reduce)
rusty1s's avatar
rusty1s committed
139
140
    elif isinstance(other, SparseTensor):
        return spspmm(src, other, reduce)
rusty1s's avatar
rusty1s committed
141
    raise ValueError
rusty1s's avatar
matmul  
rusty1s committed
142
143


rusty1s's avatar
rusty1s committed
144
145
SparseTensor.spmm = lambda self, other, reduce="sum": spmm(self, other, reduce)
SparseTensor.spspmm = lambda self, other, reduce="sum": spspmm(
rusty1s's avatar
rusty1s committed
146
    self, other, reduce)
rusty1s's avatar
rusty1s committed
147
SparseTensor.matmul = lambda self, other, reduce="sum": matmul(
rusty1s's avatar
matmul  
rusty1s committed
148
149
    self, other, reduce)
SparseTensor.__matmul__ = lambda self, other: matmul(self, other, 'sum')