aboutsummaryrefslogtreecommitdiffhomepage
path: root/tensorflow/python/kernel_tests/matrix_solve_ls_op_test.py
blob: 225a10e117fca43aad03a298556b68c3a9bd9341 (plain)
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
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
# Copyright 2015 The TensorFlow 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.
# ==============================================================================
"""Tests for tensorflow.ops.math_ops.matrix_solve."""

from __future__ import absolute_import
from __future__ import division
from __future__ import print_function

import numpy as np

from tensorflow.python.client import session
from tensorflow.python.framework import constant_op
from tensorflow.python.framework import dtypes
from tensorflow.python.framework import ops
from tensorflow.python.ops import array_ops
from tensorflow.python.ops import control_flow_ops
from tensorflow.python.ops import linalg_ops
from tensorflow.python.ops import math_ops
from tensorflow.python.ops import variables
from tensorflow.python.platform import benchmark
from tensorflow.python.platform import test as test_lib


def _AddTest(test, op_name, testcase_name, fn):
  test_name = "_".join(["test", op_name, testcase_name])
  if hasattr(test, test_name):
    raise RuntimeError("Test %s defined more than once" % test_name)
  setattr(test, test_name, fn)


def _GenerateTestData(matrix_shape, num_rhs):
  batch_shape = matrix_shape[:-2]
  matrix_shape = matrix_shape[-2:]
  m = matrix_shape[-2]
  np.random.seed(1)
  matrix = np.random.uniform(
      low=-1.0, high=1.0,
      size=np.prod(matrix_shape)).reshape(matrix_shape).astype(np.float32)
  rhs = np.ones([m, num_rhs]).astype(np.float32)
  matrix = variables.Variable(
      np.tile(matrix, batch_shape + (1, 1)), trainable=False)
  rhs = variables.Variable(np.tile(rhs, batch_shape + (1, 1)), trainable=False)
  return matrix, rhs


def _SolveWithNumpy(matrix, rhs, l2_regularizer=0):
  if l2_regularizer == 0:
    np_ans, _, _, _ = np.linalg.lstsq(matrix, rhs)
    return np_ans
  else:
    rows = matrix.shape[-2]
    cols = matrix.shape[-1]
    if rows >= cols:
      preconditioner = l2_regularizer * np.identity(cols)
      gramian = np.dot(np.conj(matrix.T), matrix) + preconditioner
      rhs = np.dot(np.conj(matrix.T), rhs)
      return np.linalg.solve(gramian, rhs)
    else:
      preconditioner = l2_regularizer * np.identity(rows)
      gramian = np.dot(matrix, np.conj(matrix.T)) + preconditioner
      z = np.linalg.solve(gramian, rhs)
      return np.dot(np.conj(matrix.T), z)


class MatrixSolveLsOpTest(test_lib.TestCase):

  def _verifySolve(self,
                   x,
                   y,
                   dtype,
                   use_placeholder,
                   fast,
                   l2_regularizer,
                   batch_shape=()):
    if not fast and l2_regularizer != 0:
      # The slow path does not support regularization.
      return
    maxdim = np.max(x.shape)
    if dtype == np.float32 or dtype == np.complex64:
      tol = maxdim * 5e-4
    else:
      tol = maxdim * 5e-7
      a = x.astype(dtype)
      b = y.astype(dtype)
      if dtype in [np.complex64, np.complex128]:
        a.imag = a.real
        b.imag = b.real
      # numpy.linalg.lstqr does not batching, so we just solve a single system
      # and replicate the solution. and residual norm.
      np_ans = _SolveWithNumpy(x, y, l2_regularizer=l2_regularizer)
      np_r = np.dot(np.conj(a.T), b - np.dot(a, np_ans))
      np_r_norm = np.sqrt(np.sum(np.conj(np_r) * np_r))
      if batch_shape is not ():
        a = np.tile(a, batch_shape + (1, 1))
        b = np.tile(b, batch_shape + (1, 1))
        np_ans = np.tile(np_ans, batch_shape + (1, 1))
        np_r_norm = np.tile(np_r_norm, batch_shape)
      with self.test_session(use_gpu=fast) as sess:
        if use_placeholder:
          a_ph = array_ops.placeholder(dtypes.as_dtype(dtype))
          b_ph = array_ops.placeholder(dtypes.as_dtype(dtype))
          feed_dict = {a_ph: a, b_ph: b}
          tf_ans = linalg_ops.matrix_solve_ls(
              a_ph, b_ph, fast=fast, l2_regularizer=l2_regularizer)
        else:
          tf_ans = linalg_ops.matrix_solve_ls(
              a, b, fast=fast, l2_regularizer=l2_regularizer)
          feed_dict = {}
          self.assertEqual(np_ans.shape, tf_ans.get_shape())
        if l2_regularizer == 0:
          # The least squares solution should satisfy A^H * (b - A*x) = 0.
          tf_r = b - math_ops.matmul(a, tf_ans)
          tf_r = math_ops.matmul(a, tf_r, adjoint_a=True)
          tf_r_norm = linalg_ops.norm(tf_r, ord="fro", axis=[-2, -1])
          tf_ans_val, tf_r_norm_val = sess.run(
              [tf_ans, tf_r_norm], feed_dict=feed_dict)
          self.assertAllClose(np_r_norm, tf_r_norm_val, atol=tol, rtol=tol)
        else:
          tf_ans_val = sess.run(tf_ans, feed_dict=feed_dict)

      self.assertEqual(np_ans.shape, tf_ans_val.shape)
      self.assertAllClose(np_ans, tf_ans_val, atol=2 * tol, rtol=2 * tol)

  def testWrongDimensions(self):
    # The matrix and right-hand sides should have the same number of rows.
    with self.test_session(use_gpu=True):
      matrix = constant_op.constant([[1., 0.], [0., 1.]])
      rhs = constant_op.constant([[1., 0.]])
      with self.assertRaises(ValueError):
        linalg_ops.matrix_solve_ls(matrix, rhs)

  def testEmpty(self):
    full = np.array([[1., 2.], [3., 4.], [5., 6.]])
    empty0 = np.empty([3, 0])
    empty1 = np.empty([0, 2])
    for fast in [True, False]:
      with self.test_session(use_gpu=True):
        tf_ans = linalg_ops.matrix_solve_ls(empty0, empty0, fast=fast).eval()
        self.assertEqual(tf_ans.shape, (0, 0))
        tf_ans = linalg_ops.matrix_solve_ls(empty0, full, fast=fast).eval()
        self.assertEqual(tf_ans.shape, (0, 2))
        tf_ans = linalg_ops.matrix_solve_ls(full, empty0, fast=fast).eval()
        self.assertEqual(tf_ans.shape, (2, 0))
        tf_ans = linalg_ops.matrix_solve_ls(empty1, empty1, fast=fast).eval()
        self.assertEqual(tf_ans.shape, (2, 2))

  def testBatchResultSize(self):
    # 3x3x3 matrices, 3x3x1 right-hand sides.
    matrix = np.array([1., 2., 3., 4., 5., 6., 7., 8., 9.] * 3).reshape(3, 3, 3)
    rhs = np.array([1., 2., 3.] * 3).reshape(3, 3, 1)
    answer = linalg_ops.matrix_solve(matrix, rhs)
    ls_answer = linalg_ops.matrix_solve_ls(matrix, rhs)
    self.assertEqual(ls_answer.get_shape(), [3, 3, 1])
    self.assertEqual(answer.get_shape(), [3, 3, 1])


def _GetSmallMatrixSolveLsOpTests(dtype, use_placeholder, fast, l2_regularizer):

  def Square(self):
    # 2x2 matrices, 2x3 right-hand sides.
    matrix = np.array([[1., 2.], [3., 4.]])
    rhs = np.array([[1., 0., 1.], [0., 1., 1.]])
    for batch_shape in (), (2, 3):
      self._verifySolve(
          matrix,
          rhs,
          dtype,
          use_placeholder,
          fast,
          l2_regularizer,
          batch_shape=batch_shape)

  def Overdetermined(self):
    # 2x2 matrices, 2x3 right-hand sides.
    matrix = np.array([[1., 2.], [3., 4.], [5., 6.]])
    rhs = np.array([[1., 0., 1.], [0., 1., 1.], [1., 1., 0.]])
    for batch_shape in (), (2, 3):
      self._verifySolve(
          matrix,
          rhs,
          dtype,
          use_placeholder,
          fast,
          l2_regularizer,
          batch_shape=batch_shape)

  def Underdetermined(self):
    # 2x2 matrices, 2x3 right-hand sides.
    matrix = np.array([[1., 2., 3], [4., 5., 6.]])
    rhs = np.array([[1., 0., 1.], [0., 1., 1.]])
    for batch_shape in (), (2, 3):
      self._verifySolve(
          matrix,
          rhs,
          dtype,
          use_placeholder,
          fast,
          l2_regularizer,
          batch_shape=batch_shape)

  return (Square, Overdetermined, Underdetermined)


def _GetLargeMatrixSolveLsOpTests(dtype, use_placeholder, fast, l2_regularizer):

  def LargeBatchSquare(self):
    np.random.seed(1)
    num_rhs = 1
    matrix_shape = (127, 127)
    matrix = np.random.uniform(
        low=-1.0, high=1.0,
        size=np.prod(matrix_shape)).reshape(matrix_shape).astype(np.float32)
    rhs = np.ones([matrix_shape[0], num_rhs]).astype(np.float32)
    self._verifySolve(
        matrix,
        rhs,
        dtype,
        use_placeholder,
        fast,
        l2_regularizer,
        batch_shape=(16, 8))

  def LargeBatchOverdetermined(self):
    np.random.seed(1)
    num_rhs = 1
    matrix_shape = (127, 64)
    matrix = np.random.uniform(
        low=-1.0, high=1.0,
        size=np.prod(matrix_shape)).reshape(matrix_shape).astype(np.float32)
    rhs = np.ones([matrix_shape[0], num_rhs]).astype(np.float32)
    self._verifySolve(
        matrix,
        rhs,
        dtype,
        use_placeholder,
        fast,
        l2_regularizer,
        batch_shape=(16, 8))

  def LargeBatchUnderdetermined(self):
    np.random.seed(1)
    num_rhs = 1
    matrix_shape = (64, 127)
    matrix = np.random.uniform(
        low=-1.0, high=1.0,
        size=np.prod(matrix_shape)).reshape(matrix_shape).astype(np.float32)
    rhs = np.ones([matrix_shape[0], num_rhs]).astype(np.float32)
    self._verifySolve(
        matrix,
        rhs,
        dtype,
        use_placeholder,
        fast,
        l2_regularizer,
        batch_shape=(16, 8))

  return (LargeBatchSquare, LargeBatchOverdetermined, LargeBatchUnderdetermined)


class MatrixSolveLsBenchmark(test_lib.Benchmark):

  matrix_shapes = [
      (4, 4),
      (8, 4),
      (4, 8),
      (10, 10),
      (10, 8),
      (8, 10),
      (16, 16),
      (16, 10),
      (10, 16),
      (101, 101),
      (101, 31),
      (31, 101),
      (256, 256),
      (256, 200),
      (200, 256),
      (1001, 1001),
      (1001, 501),
      (501, 1001),
      (1024, 1024),
      (1024, 128),
      (128, 1024),
      (2048, 2048),
      (2048, 64),
      (64, 2048),
      (513, 4, 4),
      (513, 4, 2),
      (513, 2, 4),
      (513, 16, 16),
      (513, 16, 10),
      (513, 10, 16),
      (513, 256, 256),
      (513, 256, 128),
      (513, 128, 256),
  ]

  def benchmarkMatrixSolveLsOp(self):
    run_gpu_test = test_lib.is_gpu_available(True)
    regularizer = 1.0
    for matrix_shape in self.matrix_shapes:
      for num_rhs in 1, 2, matrix_shape[-1]:

        with ops.Graph().as_default(), \
            session.Session(config=benchmark.benchmark_config()) as sess, \
            ops.device("/cpu:0"):
          matrix, rhs = _GenerateTestData(matrix_shape, num_rhs)
          x = linalg_ops.matrix_solve_ls(matrix, rhs, regularizer)
          variables.global_variables_initializer().run()
          self.run_op_benchmark(
              sess,
              control_flow_ops.group(x),
              min_iters=25,
              store_memory_usage=False,
              name=("matrix_solve_ls_cpu_shape_{matrix_shape}_num_rhs_{num_rhs}"
                   ).format(matrix_shape=matrix_shape, num_rhs=num_rhs))

        if run_gpu_test and (len(matrix_shape) < 3 or matrix_shape[0] < 513):
          with ops.Graph().as_default(), \
                session.Session(config=benchmark.benchmark_config()) as sess, \
                ops.device("/gpu:0"):
            matrix, rhs = _GenerateTestData(matrix_shape, num_rhs)
            x = linalg_ops.matrix_solve_ls(matrix, rhs, regularizer)
            variables.global_variables_initializer().run()
            self.run_op_benchmark(
                sess,
                control_flow_ops.group(x),
                min_iters=25,
                store_memory_usage=False,
                name=("matrix_solve_ls_gpu_shape_{matrix_shape}_num_rhs_"
                      "{num_rhs}").format(
                          matrix_shape=matrix_shape, num_rhs=num_rhs))


if __name__ == "__main__":
  for dtype_ in [np.float32, np.float64, np.complex64, np.complex128]:
    for use_placeholder_ in [True, False]:
      for fast_ in [True, False]:
        l2_regularizers = [0] if dtype_ == np.complex128 else [0, 0.1]
        for l2_regularizer_ in l2_regularizers:
          for test_case in _GetSmallMatrixSolveLsOpTests(
              dtype_, use_placeholder_, fast_, l2_regularizer_):
            name = "%s_%s_placeholder_%s_fast_%s_regu_%s" % (test_case.__name__,
                                                             dtype_.__name__,
                                                             use_placeholder_,
                                                             fast_,
                                                             l2_regularizer_)
            _AddTest(MatrixSolveLsOpTest, "MatrixSolveLsOpTest", name,
                     test_case)
  for dtype_ in [np.float32, np.float64, np.complex64, np.complex128]:
    for test_case in _GetLargeMatrixSolveLsOpTests(dtype_, False, True, 0.0):
      name = "%s_%s" % (test_case.__name__, dtype_.__name__)
      _AddTest(MatrixSolveLsOpTest, "MatrixSolveLsOpTest", name, test_case)

  test_lib.main()